Skip to content

Commit d5c29ff

Browse files
Cortex_m backend: Add cortex_m tester + test_add (#14510)
Note that tests are currently failing, comparing for example the call to arm_elementwise_add_s8 in op_quantized_add.cpp https://github.com/pytorch/executorch/blob/ab3100715afd21e5f0cee48675d9187152775d86/backends/cortex_m/ops/op_quantized_add.cpp#L88 with the definition in CMSIS-NN https://github.com/ARM-software/CMSIS-NN/blob/88f1982a69c00ed13dd633a63da1009c48abbb4d/Include/arm_nnfunctions.h#L1923 it seems that the args are listed in the wrong order. This will be fixed in a future patch. Minor fixes to get this to work: - Add init file to make test names unique - Update conftest to not crash is_option_enabled for tests running from external folder Signed-off-by: Adrian Lundell <[email protected]>
1 parent 98b1052 commit d5c29ff

File tree

4 files changed

+280
-1
lines changed

4 files changed

+280
-1
lines changed

backends/arm/test/conftest.py

Lines changed: 1 addition & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -118,7 +118,7 @@ def is_option_enabled(option: str, fail_if_not_enabled: bool = False) -> bool:
118118
a RuntimeError instead of returning False.
119119
"""
120120

121-
if option in pytest._test_options and pytest._test_options[option]: # type: ignore[attr-defined]
121+
if hasattr(pytest, "_test_options") and option in pytest._test_options and pytest._test_options[option]: # type: ignore[attr-defined]
122122
return True
123123
else:
124124
if fail_if_not_enabled:
Lines changed: 4 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,4 @@
1+
# Copyright 2025 Arm Limited and/or its affiliates.
2+
#
3+
# This source code is licensed under the BSD-style license found in the
4+
# LICENSE file in the root directory of this source tree.
Lines changed: 175 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,175 @@
1+
# Copyright 2025 Arm Limited and/or its affiliates.
2+
#
3+
# This source code is licensed under the BSD-style license found in the
4+
# LICENSE file in the root directory of this source tree.
5+
6+
7+
import torch
8+
from executorch.backends.arm.test.common import parametrize
9+
from executorch.backends.cortex_m.test.tester import CortexMTester, McuTestCase
10+
from executorch.backends.test.suite.operators.test_add import Model, ModelAlpha
11+
12+
13+
class CortexMSelfAdd(torch.nn.Module):
14+
ops_before_transforms = {
15+
"executorch_exir_dialects_edge__ops_aten_add_Tensor": 1,
16+
"executorch_exir_dialects_edge__ops_quantized_decomposed_quantize_per_tensor_default": 2,
17+
"executorch_exir_dialects_edge__ops_quantized_decomposed_dequantize_per_tensor_default": 2,
18+
}
19+
20+
ops_after_transforms = {
21+
"executorch_exir_dialects_edge__ops_cortex_m_quantized_add_default": 1,
22+
"executorch_exir_dialects_edge__ops_cortex_m_quantize_per_tensor_default": 1,
23+
"executorch_exir_dialects_edge__ops_cortex_m_dequantize_per_tensor_default": 1,
24+
}
25+
26+
def forward(self, x):
27+
return x + x
28+
29+
30+
class CortexMScalarAdd(Model):
31+
ops_before_transforms = {
32+
"executorch_exir_dialects_edge__ops_aten_add_Tensor": 1,
33+
"executorch_exir_dialects_edge__ops_quantized_decomposed_quantize_per_tensor_default": 2,
34+
"executorch_exir_dialects_edge__ops_quantized_decomposed_dequantize_per_tensor_default": 3,
35+
}
36+
37+
ops_after_transforms = {
38+
"executorch_exir_dialects_edge__ops_cortex_m_quantized_add_default": 1,
39+
"executorch_exir_dialects_edge__ops_cortex_m_quantize_per_tensor_default": 1,
40+
"executorch_exir_dialects_edge__ops_cortex_m_dequantize_per_tensor_default": 1,
41+
}
42+
43+
44+
class CortexMTensorAdd(Model):
45+
ops_before_transforms = {
46+
"executorch_exir_dialects_edge__ops_aten_add_Tensor": 1,
47+
"executorch_exir_dialects_edge__ops_quantized_decomposed_quantize_per_tensor_default": 3,
48+
"executorch_exir_dialects_edge__ops_quantized_decomposed_dequantize_per_tensor_default": 3,
49+
}
50+
51+
ops_after_transforms = {
52+
"executorch_exir_dialects_edge__ops_cortex_m_quantized_add_default": 1,
53+
"executorch_exir_dialects_edge__ops_cortex_m_quantize_per_tensor_default": 2,
54+
"executorch_exir_dialects_edge__ops_cortex_m_dequantize_per_tensor_default": 1,
55+
}
56+
57+
58+
class CortexMAlphaAdd(ModelAlpha):
59+
ops_before_transforms = {
60+
"executorch_exir_dialects_edge__ops_aten_add_Tensor": 1,
61+
"executorch_exir_dialects_edge__ops_quantized_decomposed_quantize_per_tensor_default": 3,
62+
"executorch_exir_dialects_edge__ops_quantized_decomposed_dequantize_per_tensor_default": 3,
63+
}
64+
65+
ops_after_transforms = {
66+
"executorch_exir_dialects_edge__ops_cortex_m_quantized_add_default": 1,
67+
"executorch_exir_dialects_edge__ops_cortex_m_quantize_per_tensor_default": 2,
68+
"executorch_exir_dialects_edge__ops_cortex_m_dequantize_per_tensor_default": 1,
69+
}
70+
71+
72+
test_cases = {
73+
"self_scalar": McuTestCase(
74+
CortexMSelfAdd(),
75+
(10.0,),
76+
),
77+
"self_rank_1": McuTestCase(
78+
CortexMSelfAdd(),
79+
(torch.linspace(-5, 5, 10),),
80+
),
81+
"self_rank_2_pos": McuTestCase(
82+
CortexMSelfAdd(),
83+
(torch.linspace(0, 1000, 10).reshape((10, 1)),),
84+
),
85+
"self_rank_3_neg": McuTestCase(
86+
CortexMSelfAdd(),
87+
(torch.linspace(-100, 0, 8).reshape((2, 2, 2)),),
88+
),
89+
"self_rank_4_small": McuTestCase(
90+
CortexMSelfAdd(),
91+
(torch.linspace(-0.1, 0.1, 16).reshape(2, 2, 2, 2),),
92+
),
93+
"self_rank_5": McuTestCase(
94+
CortexMSelfAdd(),
95+
(torch.linspace(-5, 5, 32).reshape(2, 2, 2, 2, 2),),
96+
),
97+
"scalar_scalar": McuTestCase(
98+
CortexMScalarAdd(),
99+
(-0.5, 1.0),
100+
),
101+
"tensor_scalar": McuTestCase(
102+
CortexMScalarAdd(),
103+
(torch.ones(2, 2), 1.0),
104+
),
105+
"scalar_tensor": McuTestCase(
106+
CortexMScalarAdd(),
107+
(1000.0, torch.ones(2, 2)),
108+
),
109+
"broadcast_1": McuTestCase(
110+
CortexMTensorAdd(),
111+
(torch.ones(1), torch.ones(2, 2, 2, 2)),
112+
),
113+
"broadcast_2": McuTestCase(
114+
CortexMTensorAdd(),
115+
(torch.ones((2, 1, 1, 1)), torch.ones(1)),
116+
),
117+
"broadcast_3": McuTestCase(
118+
CortexMTensorAdd(),
119+
(
120+
torch.linspace(-2, 2, 4).reshape(2, 1, 2, 1),
121+
torch.linspace(-5, 5, 4).reshape(1, 2, 1, 2),
122+
),
123+
),
124+
"alpha": McuTestCase(
125+
CortexMAlphaAdd(0.5),
126+
(
127+
torch.linspace(-10, 10, 20).reshape(4, 5),
128+
torch.linspace(-20, 20, 20).reshape(4, 5),
129+
),
130+
),
131+
}
132+
133+
134+
dialect_xfails = {
135+
"self_scalar": ("'float' object has no attribute 'fake_mode'", AttributeError),
136+
"self_rank_1": ("Output 0 does not match reference output", AssertionError),
137+
"self_rank_2_pos": ("Output 0 does not match reference output", AssertionError),
138+
"self_rank_3_neg": ("Output 0 does not match reference output", AssertionError),
139+
"self_rank_4_small": ("Output 0 does not match reference output", AssertionError),
140+
"self_rank_5": ("Output 0 does not match reference output", AssertionError),
141+
"scalar_scalar": ("'float' object has no attribute 'fake_mode'", AttributeError),
142+
"broadcast_3": ("Output 0 does not match reference output", AssertionError),
143+
"alpha": ("Expecting kwargs for aten op IR to be empty", AssertionError),
144+
}
145+
146+
147+
@parametrize("test_case", test_cases, xfails=dialect_xfails)
148+
def test_dialect_add(test_case):
149+
tester = CortexMTester(test_case.model, test_case.example_inputs)
150+
tester.test_dialect(
151+
test_case.model.ops_before_transforms, test_case.model.ops_after_transforms
152+
)
153+
154+
155+
implementation_xfails = {
156+
"self_scalar": ("'float' object has no attribute 'fake_mode'", AttributeError),
157+
"self_rank_1": ("Output 0 does not match reference output", AssertionError),
158+
"self_rank_2_pos": ("Output 0 does not match reference output", AssertionError),
159+
"self_rank_3_neg": ("Output 0 does not match reference output", AssertionError),
160+
"self_rank_4_small": ("Output 0 does not match reference output", AssertionError),
161+
"self_rank_5": ("Output 0 does not match reference output", AssertionError),
162+
"scalar_scalar": ("'float' object has no attribute 'fake_mode'", AttributeError),
163+
"tensor_scalar": ("Output 0 does not match reference output", AssertionError),
164+
"scalar_tensor": ("Output 0 does not match reference output", AssertionError),
165+
"broadcast_1": ("Output 0 does not match reference output", AssertionError),
166+
"broadcast_2": ("Output 0 does not match reference output", AssertionError),
167+
"broadcast_3": ("Output 0 does not match reference output", AssertionError),
168+
"alpha": ("Expecting kwargs for aten op IR to be empty", AssertionError),
169+
}
170+
171+
172+
@parametrize("test_case", test_cases, xfails=implementation_xfails)
173+
def test_implementation_add(test_case):
174+
tester = CortexMTester(test_case.model, test_case.example_inputs)
175+
tester.test_implementation()

backends/cortex_m/test/tester.py

Lines changed: 100 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,100 @@
1+
# Copyright 2025 Arm Limited and/or its affiliates.
2+
#
3+
# This source code is licensed under the BSD-style license found in the
4+
# LICENSE file in the root directory of this source tree.
5+
6+
7+
from dataclasses import dataclass
8+
from typing import Any
9+
10+
import torch
11+
12+
from backends.xnnpack.quantizer.xnnpack_quantizer import (
13+
get_symmetric_quantization_config,
14+
XNNPACKQuantizer,
15+
)
16+
from executorch.backends.arm.test.common import get_u55_compile_spec
17+
from executorch.backends.arm.test.tester.arm_tester import Serialize
18+
from executorch.backends.cortex_m.passes.quantized_op_fusion_pass import (
19+
QuantizedOpFusionPass,
20+
)
21+
22+
from executorch.backends.cortex_m.passes.replace_quant_nodes_pass import (
23+
ReplaceQuantNodesPass,
24+
)
25+
from executorch.backends.test.harness import Tester as TesterBase
26+
from executorch.backends.test.harness.stages import (
27+
Export,
28+
Quantize,
29+
RunPasses,
30+
StageType,
31+
ToEdgeTransformAndLower,
32+
ToExecutorch,
33+
)
34+
from executorch.backends.xnnpack._passes import XNNPACKPassManager
35+
36+
37+
class CortexMQuantize(Quantize):
38+
def __init__(self):
39+
quantizer = XNNPACKQuantizer()
40+
config = get_symmetric_quantization_config()
41+
super().__init__(quantizer, config)
42+
43+
44+
class CortexMRunPasses(RunPasses):
45+
def __init__(self):
46+
super().__init__(
47+
XNNPACKPassManager, pass_list=[QuantizedOpFusionPass, ReplaceQuantNodesPass]
48+
)
49+
50+
51+
class CortexMSerialize(Serialize):
52+
def __init__(self):
53+
compile_spec = get_u55_compile_spec()
54+
super().__init__(compile_spec, 1024)
55+
56+
57+
cortex_m_stage_classes = {
58+
StageType.EXPORT: Export,
59+
StageType.QUANTIZE: CortexMQuantize,
60+
StageType.RUN_PASSES: CortexMRunPasses,
61+
StageType.SERIALIZE: Serialize,
62+
StageType.TO_EDGE_TRANSFORM_AND_LOWER: ToEdgeTransformAndLower,
63+
StageType.TO_EXECUTORCH: ToExecutorch,
64+
StageType.SERIALIZE: CortexMSerialize,
65+
}
66+
67+
68+
class CortexMTester(TesterBase):
69+
def __init__(self, module, example_inputs):
70+
super().__init__(module, example_inputs, cortex_m_stage_classes)
71+
72+
def test_dialect(self, ops_before_transforms, ops_after_transforms, qtol=0):
73+
"""
74+
Test the python dialect op implementation.
75+
"""
76+
self.quantize()
77+
self.export()
78+
self.to_edge_transform_and_lower()
79+
self.check_count(ops_before_transforms)
80+
self.run_passes()
81+
self.check_count(ops_after_transforms)
82+
self.run_method_and_compare_outputs(inputs=self.example_inputs, qtol=qtol)
83+
84+
def test_implementation(self, qtol=0):
85+
"""
86+
Test the optimized op implementation in simulation
87+
"""
88+
self.quantize()
89+
self.export()
90+
self.to_edge_transform_and_lower()
91+
self.run_passes()
92+
self.to_executorch()
93+
self.serialize()
94+
self.run_method_and_compare_outputs(inputs=self.example_inputs, qtol=qtol)
95+
96+
97+
@dataclass
98+
class McuTestCase:
99+
model: torch.nn.Module
100+
example_inputs: tuple[Any]

0 commit comments

Comments
 (0)