Skip to content

Commit

Permalink
Fix ONNX operator_export_type on the new registry
Browse files Browse the repository at this point in the history
  • Loading branch information
thiagocrepaldi committed Oct 26, 2022
1 parent 44d7ba7 commit 06dc424
Show file tree
Hide file tree
Showing 2 changed files with 112 additions and 14 deletions.
85 changes: 85 additions & 0 deletions test/onnx/test_pytorch_onnx_no_runtime.py
Expand Up @@ -777,6 +777,91 @@ def forward(self, x):
model, inputs, f, dynamic_axes={"x": [0, 1]}, input_names=["x"]
)

@common_utils.skipIfNoCaffe2
def test_caffe2_aten_fallback(self):
class ModelWithAtenNotONNXOp(torch.nn.Module):
def forward(self, x, y):
abcd = x + y
defg = torch.linalg.qr(abcd)
return defg

x = torch.rand(3, 4)
y = torch.rand(3, 4)
f = io.BytesIO()
torch.onnx.export(
ModelWithAtenNotONNXOp(),
(x, y),
f,
do_constant_folding=False,
operator_export_type=torch.onnx.OperatorExportTypes.ONNX_ATEN_FALLBACK,
# support for linalg.qr was added in later op set versions.
opset_version=9,
)
onnx_model = onnx.load(io.BytesIO(f.getvalue()))
self.assertAtenOp(onnx_model, "linalg_qr")

@common_utils.skipIfNoCaffe2
def test_caffe2_onnx_aten(self):
class ModelWithAtenFmod(torch.nn.Module):
def forward(self, x, y):
return torch.fmod(x, y)

x = torch.randn(3, 4, dtype=torch.float32)
y = torch.randn(3, 4, dtype=torch.float32)
f = io.BytesIO()
torch.onnx.export(
ModelWithAtenFmod(),
(x, y),
f,
do_constant_folding=False,
operator_export_type=torch.onnx.OperatorExportTypes.ONNX_ATEN,
opset_version=10, # or higher
)
onnx_model = onnx.load(io.BytesIO(f.getvalue()))
assert onnx_model.graph.node[0].op_type == "Mod"

@common_utils.skipIfCaffe2
def test_aten_fallback(self):
class ModelWithAtenNotONNXOp(torch.nn.Module):
def forward(self, x, y):
abcd = x + y
defg = torch.linalg.qr(abcd)
return defg

x = torch.rand(3, 4)
y = torch.rand(3, 4)
f = io.BytesIO()
torch.onnx.export(
ModelWithAtenNotONNXOp(),
(x, y),
f,
do_constant_folding=False,
operator_export_type=torch.onnx.OperatorExportTypes.ONNX_ATEN_FALLBACK,
# support for linalg.qr was added in later op set versions.
opset_version=9,
)
onnx_model = onnx.load(io.BytesIO(f.getvalue()))
self.assertAtenOp(onnx_model, "linalg_qr")

@common_utils.skipIfCaffe2
def test_onnx_aten(self):
class ModelWithAtenFmod(torch.nn.Module):
def forward(self, x, y):
return torch.fmod(x, y)

x = torch.randn(3, 4, dtype=torch.float32)
y = torch.randn(3, 4, dtype=torch.float32)
f = io.BytesIO()
torch.onnx.export(
ModelWithAtenFmod(),
(x, y),
f,
do_constant_folding=False,
operator_export_type=torch.onnx.OperatorExportTypes.ONNX_ATEN,
)
onnx_model = onnx.load(io.BytesIO(f.getvalue()))
self.assertAtenOp(onnx_model, "fmod", "Tensor")


if __name__ == "__main__":
common_utils.run_tests()
41 changes: 27 additions & 14 deletions torch/onnx/utils.py
Expand Up @@ -1748,8 +1748,8 @@ def _should_aten_fallback(
is_aten_fallback_export = (
operator_export_type == _C_onnx.OperatorExportTypes.ONNX_ATEN_FALLBACK
)
return is_onnx_aten_export or (
not is_exportable_aten_op and is_aten_fallback_export
return name.startswith("aten::") and (
is_onnx_aten_export or (not is_exportable_aten_op and is_aten_fallback_export)
)


Expand Down Expand Up @@ -1844,6 +1844,30 @@ def _run_symbolic_function(
env=env,
)

# Direct ATen export requested
# For BUILD_CAFFE2=0 builds, if domain=="aten" and operator_export_type==ONNX_ATEN,
# an aten::ATen operator is created regardless of symbolics existence
# For BUILD_CAFFE2=1, the same applies only if there is no symbolic available
if (
_C_onnx._CAFFE2_ATEN_FALLBACK
and not registration.registry.is_registered_op(ns_op_name, opset_version)
) or (
not _C_onnx._CAFFE2_ATEN_FALLBACK
and _should_aten_fallback(ns_op_name, opset_version, operator_export_type)
):
attrs = {
k + "_" + node.kindOf(k)[0]: symbolic_helper._node_get(node, k)
for k in node.attributeNames()
}
outputs = node.outputsSize()
attrs["outputs"] = outputs
return graph_context.at(
op_name,
*inputs,
overload_name=_get_aten_op_overload_name(node),
**attrs,
)

try:
# Caffe2-specific: Quantized op symbolics are registered for opset 9 only.
if symbolic_helper.is_caffe2_aten_fallback() and opset_version == 9:
Expand All @@ -1861,6 +1885,7 @@ def _run_symbolic_function(
if symbolic_function_group is not None:
symbolic_fn = symbolic_function_group.get(opset_version)
if symbolic_fn is not None:
# TODO Wrap almost identical attrs assignment or comment the difference.
attrs = {
k: symbolic_helper._node_get(node, k) for k in node.attributeNames()
}
Expand All @@ -1874,18 +1899,6 @@ def _run_symbolic_function(
# Clone node to trigger ONNX shape inference
return graph_context.op(op_name, *inputs, **attrs, outputs=node.outputsSize()) # type: ignore[attr-defined]

if _should_aten_fallback(ns_op_name, opset_version, operator_export_type):
# Direct ATen export requested
outputs = node.outputsSize()
attrs["outputs"] = outputs
# `overload_name` is set for non-Caffe2 builds only
return graph_context.at(
op_name,
*inputs,
overload_name=_get_aten_op_overload_name(node),
**attrs,
)

raise errors.UnsupportedOperatorError(
domain,
op_name,
Expand Down

0 comments on commit 06dc424

Please sign in to comment.