Skip to content

Commit d509ee3

Browse files
authored
Arm backend: Refactor models to allow for TOSA 1.0 (#10904)
### Summary Update model unit tests to use the new test infrastructure pipeline.
1 parent 379129d commit d509ee3

File tree

5 files changed

+360
-416
lines changed

5 files changed

+360
-416
lines changed

backends/arm/test/models/test_conformer.py

Lines changed: 98 additions & 102 deletions
Original file line numberDiff line numberDiff line change
@@ -3,32 +3,37 @@
33
# This source code is licensed under the BSD-style license found in the
44
# LICENSE file in the root directory of this source tree.
55

6-
import unittest
6+
from typing import Tuple
7+
8+
import pytest
79

810
import torch
9-
from executorch.backends.arm.test import common, conftest
1011

11-
from executorch.backends.arm.test.tester.arm_tester import ArmTester
12+
from executorch.backends.arm.test import common
13+
from executorch.backends.arm.test.tester.test_pipeline import (
14+
EthosU55PipelineBI,
15+
EthosU85PipelineBI,
16+
TosaPipelineBI,
17+
TosaPipelineMI,
18+
)
1219

1320
from torchaudio.models import Conformer
1421

22+
input_t = Tuple[torch.Tensor, torch.IntTensor] # Input x, y
23+
1524

1625
def get_test_inputs(dim, lengths, num_examples):
1726
return (torch.rand(num_examples, int(lengths.max()), dim), lengths)
1827

1928

20-
class TestConformer(unittest.TestCase):
29+
class TestConformer:
2130
"""Tests Torchaudio Conformer"""
2231

2332
# Adjust nbr below as we increase op support. Note: most of the delegates
2433
# calls are directly consecutive to each other in the .pte. The reason
2534
# for that is some assert ops are removed by passes in the
2635
# .to_executorch step, i.e. after Arm partitioner.
27-
ops_after_partitioner = {
28-
"executorch_exir_dialects_edge__ops_aten_max_default": 1,
29-
"torch.ops.aten._assert_scalar.default": 7,
30-
"torch.ops.aten._local_scalar_dense.default": 1,
31-
}
36+
aten_ops = ["torch.ops.aten._assert_scalar.default"]
3237

3338
dim = 16
3439
num_examples = 10
@@ -43,96 +48,87 @@ class TestConformer(unittest.TestCase):
4348
)
4449
conformer = conformer.eval()
4550

46-
def test_conformer_tosa_MI(self):
47-
(
48-
ArmTester(
49-
self.conformer,
50-
example_inputs=self.model_example_inputs,
51-
compile_spec=common.get_tosa_compile_spec(tosa_spec="TOSA-0.80+MI"),
52-
)
53-
.export()
54-
.to_edge_transform_and_lower()
55-
.dump_operator_distribution()
56-
.check_count(self.ops_after_partitioner)
57-
.to_executorch()
58-
# TODO(MLETORCH-632): Fix numerical errors
59-
.run_method_and_compare_outputs(
60-
rtol=1.0,
61-
atol=5.0,
62-
inputs=get_test_inputs(self.dim, self.lengths, self.num_examples),
63-
)
64-
)
65-
66-
def test_conformer_tosa_BI(self):
67-
(
68-
ArmTester(
69-
self.conformer,
70-
example_inputs=self.model_example_inputs,
71-
compile_spec=common.get_tosa_compile_spec(tosa_spec="TOSA-0.80+BI"),
72-
)
73-
.quantize()
74-
.export()
75-
.to_edge_transform_and_lower()
76-
.to_executorch()
77-
.run_method_and_compare_outputs(
78-
qtol=1.0,
79-
rtol=1.0,
80-
atol=5.0,
81-
inputs=get_test_inputs(self.dim, self.lengths, self.num_examples),
82-
)
83-
)
84-
85-
def test_conformer_u55_BI(self):
86-
tester = (
87-
ArmTester(
88-
self.conformer,
89-
example_inputs=self.model_example_inputs,
90-
compile_spec=common.get_u55_compile_spec(),
91-
)
92-
.quantize()
93-
.export()
94-
.to_edge_transform_and_lower()
95-
.to_executorch()
96-
.serialize()
97-
)
98-
99-
if conftest.is_option_enabled("corstone_fvp"):
100-
try:
101-
tester.run_method_and_compare_outputs(
102-
qtol=1.0,
103-
rtol=1.0,
104-
atol=5.0,
105-
inputs=get_test_inputs(self.dim, self.lengths, self.num_examples),
106-
)
107-
self.fail(
108-
"TODO(MLETORCH-635): Expected failure under FVP option, but test passed."
109-
)
110-
except Exception:
111-
pass
112-
113-
def test_conformer_u85_BI(self):
114-
tester = (
115-
ArmTester(
116-
self.conformer,
117-
example_inputs=self.model_example_inputs,
118-
compile_spec=common.get_u85_compile_spec(),
119-
)
120-
.quantize()
121-
.export()
122-
.to_edge_transform_and_lower()
123-
.to_executorch()
124-
.serialize()
125-
)
126-
if conftest.is_option_enabled("corstone_fvp"):
127-
try:
128-
tester.run_method_and_compare_outputs(
129-
qtol=1.0,
130-
rtol=1.0,
131-
atol=5.0,
132-
inputs=get_test_inputs(self.dim, self.lengths, self.num_examples),
133-
)
134-
self.fail(
135-
"TODO(MLETORCH-635): Expected failure under FVP option, but test passed."
136-
)
137-
except Exception:
138-
pass
51+
52+
def test_conformer_tosa_MI():
53+
pipeline = TosaPipelineMI[input_t](
54+
TestConformer.conformer,
55+
TestConformer.model_example_inputs,
56+
aten_op=TestConformer.aten_ops,
57+
exir_op=[],
58+
use_to_edge_transform_and_lower=True,
59+
)
60+
pipeline.change_args(
61+
"run_method_and_compare_outputs",
62+
get_test_inputs(
63+
TestConformer.dim, TestConformer.lengths, TestConformer.num_examples
64+
),
65+
rtol=1.0,
66+
atol=5.0,
67+
)
68+
pipeline.run()
69+
70+
71+
def test_conformer_tosa_BI():
72+
pipeline = TosaPipelineBI[input_t](
73+
TestConformer.conformer,
74+
TestConformer.model_example_inputs,
75+
aten_op=TestConformer.aten_ops,
76+
exir_op=[],
77+
use_to_edge_transform_and_lower=True,
78+
)
79+
pipeline.pop_stage("check_count.exir")
80+
pipeline.change_args(
81+
"run_method_and_compare_outputs",
82+
get_test_inputs(
83+
TestConformer.dim, TestConformer.lengths, TestConformer.num_examples
84+
),
85+
rtol=1.0,
86+
atol=5.0,
87+
)
88+
pipeline.run()
89+
90+
91+
@common.XfailIfNoCorstone300
92+
@pytest.mark.xfail(
93+
reason="TODO(MLETORCH-635): Expected failure under FVP option, but test passed."
94+
)
95+
def test_conformer_u55_BI():
96+
pipeline = EthosU55PipelineBI[input_t](
97+
TestConformer.conformer,
98+
TestConformer.model_example_inputs,
99+
aten_ops=TestConformer.aten_ops,
100+
exir_ops=[],
101+
use_to_edge_transform_and_lower=True,
102+
run_on_fvp=True,
103+
)
104+
pipeline.change_args(
105+
"run_method_and_compare_outputs",
106+
get_test_inputs(
107+
TestConformer.dim, TestConformer.lengths, TestConformer.num_examples
108+
),
109+
rtol=1.0,
110+
atol=5.0,
111+
)
112+
pipeline.run()
113+
114+
115+
@common.XfailIfNoCorstone320
116+
@pytest.mark.xfail(reason="All IO needs to have the same data type (MLETORCH-635)")
117+
def test_conformer_u85_BI():
118+
pipeline = EthosU85PipelineBI[input_t](
119+
TestConformer.conformer,
120+
TestConformer.model_example_inputs,
121+
aten_ops=TestConformer.aten_ops,
122+
exir_ops=[],
123+
use_to_edge_transform_and_lower=True,
124+
run_on_fvp=True,
125+
)
126+
pipeline.change_args(
127+
"run_method_and_compare_outputs",
128+
get_test_inputs(
129+
TestConformer.dim, TestConformer.lengths, TestConformer.num_examples
130+
),
131+
rtol=1.0,
132+
atol=5.0,
133+
)
134+
pipeline.run()

backends/arm/test/models/test_dl3_arm.py

Lines changed: 71 additions & 76 deletions
Original file line numberDiff line numberDiff line change
@@ -3,92 +3,87 @@
33
# This source code is licensed under the BSD-style license found in the
44
# LICENSE file in the root directory of this source tree.
55

6-
import unittest
6+
from typing import Tuple
77

88
import pytest
99

10-
from executorch.backends.arm.test import common, conftest
10+
import torch
11+
12+
from executorch.backends.arm.test import common
13+
14+
from executorch.backends.arm.test.tester.test_pipeline import (
15+
EthosU55PipelineBI,
16+
EthosU85PipelineBI,
17+
TosaPipelineBI,
18+
TosaPipelineMI,
19+
)
1120

12-
from executorch.backends.arm.test.tester.arm_tester import ArmTester
1321
from executorch.examples.models import deeplab_v3
1422

23+
input_t = Tuple[torch.Tensor] # Input x
24+
1525

16-
class TestDl3(unittest.TestCase):
26+
class TestDl3:
1727
"""Tests DeepLabv3."""
1828

1929
dl3 = deeplab_v3.DeepLabV3ResNet50Model()
2030
model_example_inputs = dl3.get_example_inputs()
2131
dl3 = dl3.get_eager_model()
2232

23-
@unittest.expectedFailure
24-
def test_dl3_tosa_MI(self):
25-
(
26-
ArmTester(
27-
self.dl3,
28-
example_inputs=self.model_example_inputs,
29-
compile_spec=common.get_tosa_compile_spec("TOSA-0.80+MI"),
30-
)
31-
.export()
32-
.to_edge_transform_and_lower()
33-
.to_executorch()
34-
.run_method_and_compare_outputs(inputs=self.dl3.get_example_inputs())
35-
)
36-
37-
@unittest.expectedFailure
38-
def test_dl3_tosa_BI(self):
39-
(
40-
ArmTester(
41-
self.dl3,
42-
example_inputs=self.model_example_inputs,
43-
compile_spec=common.get_tosa_compile_spec("TOSA-0.80+BI"),
44-
)
45-
.quantize()
46-
.export()
47-
.to_edge_transform_and_lower()
48-
.to_executorch()
49-
.run_method_and_compare_outputs(
50-
atol=1.0, qtol=1, inputs=self.dl3.get_example_inputs()
51-
)
52-
)
53-
54-
@pytest.mark.slow
55-
@pytest.mark.corstone_fvp
56-
@unittest.skip
57-
def test_dl3_u55_BI(self):
58-
tester = (
59-
ArmTester(
60-
self.dl3,
61-
example_inputs=self.model_example_inputs,
62-
compile_spec=common.get_u55_compile_spec(),
63-
)
64-
.quantize()
65-
.export()
66-
.to_edge_transform_and_lower()
67-
.to_executorch()
68-
.serialize()
69-
)
70-
if conftest.is_option_enabled("corstone_fvp"):
71-
tester.run_method_and_compare_outputs(
72-
atol=1.0, qtol=1, inputs=self.dl3.get_example_inputs()
73-
)
74-
75-
@pytest.mark.slow
76-
@pytest.mark.corstone_fvp
77-
@unittest.skip
78-
def test_dl3_u85_BI(self):
79-
tester = (
80-
ArmTester(
81-
self.dl3,
82-
example_inputs=self.model_example_inputs,
83-
compile_spec=common.get_u85_compile_spec(),
84-
)
85-
.quantize()
86-
.export()
87-
.to_edge_transform_and_lower()
88-
.to_executorch()
89-
.serialize()
90-
)
91-
if conftest.is_option_enabled("corstone_fvp"):
92-
tester.run_method_and_compare_outputs(
93-
atol=1.0, qtol=1, inputs=self.dl3.get_example_inputs()
94-
)
33+
34+
def test_dl3_tosa_MI():
35+
pipeline = TosaPipelineMI[input_t](
36+
TestDl3.dl3,
37+
TestDl3.model_example_inputs,
38+
aten_op=[],
39+
exir_op=[],
40+
)
41+
pipeline.change_args(
42+
"run_method_and_compare_outputs", rtol=1.0, atol=1.0
43+
) # TODO: MLETORCH-1036 decrease tolerance
44+
pipeline.run()
45+
46+
47+
def test_dl3_tosa_BI():
48+
pipeline = TosaPipelineBI[input_t](
49+
TestDl3.dl3,
50+
TestDl3.model_example_inputs,
51+
aten_op=[],
52+
exir_op=[],
53+
)
54+
pipeline.change_args(
55+
"run_method_and_compare_outputs", rtol=1.0, atol=1.0
56+
) # TODO: MLETORCH-1036 decrease tolerance
57+
pipeline.run()
58+
59+
60+
@common.XfailIfNoCorstone300
61+
@pytest.mark.skip(reason="upsample_bilinear2d operator is not supported on U55")
62+
def test_dl3_u55_BI():
63+
pipeline = EthosU55PipelineBI[input_t](
64+
TestDl3.dl3,
65+
TestDl3.model_example_inputs,
66+
aten_ops=[],
67+
exir_ops=[],
68+
run_on_fvp=True,
69+
)
70+
pipeline.change_args(
71+
"run_method_and_compare_outputs", rtol=1.0, atol=1.0
72+
) # TODO: MLETORCH-1036 decrease tolerance
73+
pipeline.run()
74+
75+
76+
@common.XfailIfNoCorstone320
77+
@pytest.mark.skip(reason="Runs out of memory on U85")
78+
def test_dl3_u85_BI():
79+
pipeline = EthosU85PipelineBI[input_t](
80+
TestDl3.dl3,
81+
TestDl3.model_example_inputs,
82+
aten_ops=[],
83+
exir_ops=[],
84+
run_on_fvp=True,
85+
)
86+
pipeline.change_args(
87+
"run_method_and_compare_outputs", rtol=1.0, atol=1.0
88+
) # TODO: MLETORCH-1036 decrease tolerance
89+
pipeline.run()

0 commit comments

Comments
 (0)