Skip to content

Commit c404281

Browse files
committed
Convert torch._C._linalg.linalg_vector_norm to torch.linalg.vector_norm
1 parent c65f7fa commit c404281

File tree

2 files changed

+22
-2
lines changed

2 files changed

+22
-2
lines changed

graph_net/torch/backend/unstable_to_stable_backend.py

Lines changed: 21 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -147,7 +147,27 @@ def _impl_unstable_to_stable_special_logit(self, gm):
147147

148148
return gm
149149

150-
# replace this line with modification code for task 116 (torch._C._linalg.linalg_vector_norm)
150+
def _impl_unstable_to_stable_linalg_vector_norm(self, gm):
151+
"""
152+
Convert torch._C._linalg.linalg_vector_norm to torch.linalg.vector_norm
153+
"""
154+
# Update graph nodes: replace torch._C._linalg.linalg_vector_norm with torch.linalg.vector_norm
155+
issue_nodes = (
156+
node
157+
for node in gm.graph.nodes
158+
if node.op == "call_function"
159+
if hasattr(node.target, "__module__")
160+
if node.target.__module__ == "torch._C._linalg"
161+
if hasattr(node.target, "__name__")
162+
if node.target.__name__ == "linalg_vector_norm"
163+
)
164+
for node in issue_nodes:
165+
node.target = torch.linalg.vector_norm
166+
167+
# Recompile the graph
168+
gm.recompile()
169+
170+
return gm
151171

152172
# replace this line with modification code for task 117 (torch._C._linalg.linalg_norm)
153173

graph_net/torch/fx_graph_serialize_util.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -25,7 +25,7 @@ def serialize_graph_module_to_str(gm: torch.fx.GraphModule) -> str:
2525
(r"torch\._C\._fft\.fft_rfft\(", "torch.fft.rfft("),
2626
(r"torch\._C\._fft\.fft_fftn\(", "torch.fft.fftn("),
2727
(r"torch\._C\._special\.special_logit\(", "torch.special.logit("),
28-
# replace this line with modification code for task 116 (torch._C._linalg.linalg_vector_norm)
28+
(r"torch\._C\._linalg\.linalg_vector_norm\(", "torch.linalg.vector_norm("),
2929
# replace this line with modification code for task 117 (torch._C._linalg.linalg_norm)
3030
# replace this line with modification code for task 118 (torch._C._nn.softplus)
3131
# replace this line with modification code for task 119 (torch._C._nn.one_hot)

0 commit comments

Comments
 (0)