Skip to content

Commit cd343ee

Browse files
cccclaifacebook-github-bot
authored andcommitted
support qnn mean (dim=None) (pytorch#14675)
Summary: Address mean op lower failure. When dim is not specified, it will take mean across all axes. For QNN, we need to get axes based on input shape Differential Revision: D83520776
1 parent 19be2a3 commit cd343ee

File tree

3 files changed

+135
-24
lines changed

3 files changed

+135
-24
lines changed

backends/qualcomm/builders/op_mean_dim.py

Lines changed: 14 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -40,7 +40,20 @@ def define_node(
4040
)
4141

4242
# mean dims and keep dims
43-
mean_dims = cast(List[int], node.args[1])
43+
rank = len(input_node.meta["val"].shape)
44+
45+
if rank == 0:
46+
raise RuntimeError("Mean doesn't support 0d input, please report a bug in https://github.com/pytorch/executorch/issues")
47+
48+
dim_arg = node.args[1]
49+
50+
if dim_arg is None or len(dim_arg) == 0:
51+
mean_dims = list(range(rank)) # reduce over all dims
52+
elif isinstance(dim_arg, int):
53+
mean_dims = [dim_arg]
54+
else:
55+
mean_dims = list(dim_arg)
56+
print("mean_dims: ", mean_dims, "rank: ", rank)
4457
mean_dims = [
4558
mean_dim % len(input_node.meta["val"].shape) for mean_dim in mean_dims
4659
]

backends/qualcomm/tests/models.py

Lines changed: 12 additions & 12 deletions
Original file line numberDiff line numberDiff line change
@@ -5,7 +5,7 @@
55
# LICENSE file in the root directory of this source tree.
66

77
import torch
8-
8+
from typing import Optional, Union, Tuple, List
99

1010
# module with related operator only
1111

@@ -1332,20 +1332,20 @@ def forward(self, x):
13321332
return self.max_pool2d(x)
13331333

13341334

1335-
class MeanWKeppDim(torch.nn.Module):
1336-
def __init__(self):
1337-
super().__init__()
1338-
1339-
def forward(self, x):
1340-
return torch.mean(x, (-1, -2), keepdim=True)
1341-
1342-
1343-
class MeanWOKeppDim(torch.nn.Module):
1344-
def __init__(self):
1335+
class Mean(torch.nn.Module):
1336+
def __init__(
1337+
self,
1338+
dim: Optional[Union[int, Tuple[int, ...], List[int]]] = None,
1339+
keepdim: bool = False,
1340+
dtype: Optional[torch.dtype] = None,
1341+
):
13451342
super().__init__()
1343+
self.dim = dim
1344+
self.keepdim = keepdim
1345+
self.dtype = dtype
13461346

13471347
def forward(self, x):
1348-
return torch.mean(x, (-1, -2))
1348+
return torch.mean(x, dim=self.dim, keepdim=self.keepdim, dtype=self.dtype)
13491349

13501350

13511351
class MaskedFill(torch.nn.Module):

backends/qualcomm/tests/test_qnn_delegate.py

Lines changed: 109 additions & 11 deletions
Original file line numberDiff line numberDiff line change
@@ -1018,12 +1018,60 @@ def test_qnn_backend_max_pool2d(self):
10181018
sample_input = (torch.randn(4, 3, 24, 24),)
10191019
self.lower_module_and_test_output(module, sample_input)
10201020

1021-
def test_qnn_backend_mean_dim(self):
1022-
modules = [MeanWKeppDim(), MeanWOKeppDim()] # noqa: F405
1023-
sample_input = (torch.randn([2, 5, 1, 3]),)
1024-
for i, module in enumerate(modules):
1021+
def test_qnn_backend_mean(self):
1022+
test_comb = [
1023+
# Reduce over last two dims, keepdim=True
1024+
{
1025+
QCOM_MODULE: Mean(dim=(-1, -2), keepdim=True),
1026+
QCOM_SAMPLE_INPUTS: (torch.randn([2, 5, 1, 3]),),
1027+
},
1028+
# Reduce over last two dims, keepdim=False
1029+
{
1030+
QCOM_MODULE: Mean(dim=(-1, -2), keepdim=False),
1031+
QCOM_SAMPLE_INPUTS: (torch.randn([2, 5, 1, 3]),),
1032+
},
1033+
# Default: reduce all dims
1034+
{
1035+
QCOM_MODULE: Mean(),
1036+
QCOM_SAMPLE_INPUTS: (torch.randn(10, 10),),
1037+
},
1038+
# TODO: To be enabled via reshape input to 1d tensor
1039+
# # Scalar case
1040+
# {
1041+
# QCOM_MODULE: Mean(),
1042+
# QCOM_SAMPLE_INPUTS: (torch.tensor(5.0),),
1043+
# },
1044+
# Edge case: dim is a empty list
1045+
{
1046+
QCOM_MODULE: Mean(dim=[]),
1047+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
1048+
},
1049+
# Edge case: reduce along dim=0 (batch dimension)
1050+
{
1051+
QCOM_MODULE: Mean(dim=0),
1052+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
1053+
},
1054+
# Edge case: reduce along dim=0 with keepdim=True
1055+
{
1056+
QCOM_MODULE: Mean(dim=0, keepdim=True),
1057+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
1058+
},
1059+
# Edge case: reduce along multiple dims
1060+
{
1061+
QCOM_MODULE: Mean(dim=(0, 2)),
1062+
QCOM_SAMPLE_INPUTS: (torch.randn(3, 4, 5),),
1063+
},
1064+
# Edge case: high-dimensional tensor
1065+
{
1066+
QCOM_MODULE: Mean(dim=(1, 3), keepdim=True),
1067+
QCOM_SAMPLE_INPUTS: (torch.randn(2, 3, 4, 5, 6),),
1068+
},
1069+
]
1070+
1071+
for i, test in enumerate(test_comb):
10251072
with self.subTest(i=i):
1026-
self.lower_module_and_test_output(module, sample_input)
1073+
print("running test i: ", i)
1074+
self.lower_module_and_test_output(test[QCOM_MODULE], test[QCOM_SAMPLE_INPUTS])
10271075

10281076
@unittest.skip("failed to lower in QNN 2.26")
10291077
def test_qnn_backend_mha(self):
@@ -2666,13 +2714,63 @@ def test_qnn_backend_max_pool2d(self):
26662714
module = self.get_qdq_module(module, sample_input)
26672715
self.lower_module_and_test_output(module, sample_input)
26682716

2669-
def test_qnn_backend_mean_dim(self):
2670-
modules = [MeanWKeppDim(), MeanWOKeppDim()] # noqa: F405
2671-
sample_input = (torch.randn([2, 5, 1, 3]),)
2672-
for i, module in enumerate(modules):
2717+
def test_qnn_backend_mean(self):
2718+
test_comb = [
2719+
# Reduce over last two dims, keepdim=True
2720+
{
2721+
QCOM_MODULE: Mean(dim=(-1, -2), keepdim=True),
2722+
QCOM_SAMPLE_INPUTS: (torch.randn([2, 5, 1, 3]),),
2723+
},
2724+
# Reduce over last two dims, keepdim=False
2725+
{
2726+
QCOM_MODULE: Mean(dim=(-1, -2), keepdim=False),
2727+
QCOM_SAMPLE_INPUTS: (torch.randn([2, 5, 1, 3]),),
2728+
},
2729+
# Default: reduce all dims
2730+
{
2731+
QCOM_MODULE: Mean(),
2732+
QCOM_SAMPLE_INPUTS: (torch.randn(10, 10),),
2733+
},
2734+
# TODO: To be enabled via reshape input to 1d tensor
2735+
# Scalar case
2736+
# {
2737+
# QCOM_MODULE: Mean(),
2738+
# QCOM_SAMPLE_INPUTS: (torch.tensor(5.0),),
2739+
# },
2740+
# Edge case: dim is a empty list
2741+
{
2742+
QCOM_MODULE: Mean(dim=[]),
2743+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
2744+
},
2745+
# Edge case: reduce along dim=0 (batch dimension)
2746+
{
2747+
QCOM_MODULE: Mean(dim=0),
2748+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
2749+
},
2750+
# Edge case: reduce along dim=0 with keepdim=True
2751+
{
2752+
QCOM_MODULE: Mean(dim=0, keepdim=True),
2753+
QCOM_SAMPLE_INPUTS: (torch.randn(4, 6, 8),),
2754+
},
2755+
# Edge case: reduce along multiple dims
2756+
{
2757+
QCOM_MODULE: Mean(dim=(0, 2)),
2758+
QCOM_SAMPLE_INPUTS: (torch.randn(3, 4, 5),),
2759+
},
2760+
# Edge case: high-dimensional tensor
2761+
{
2762+
QCOM_MODULE: Mean(dim=(1, 3), keepdim=True),
2763+
QCOM_SAMPLE_INPUTS: (torch.randn(2, 3, 4, 5, 6),),
2764+
},
2765+
]
2766+
2767+
for i, test in enumerate(test_comb):
26732768
with self.subTest(i=i):
2674-
module = self.get_qdq_module(module, sample_input)
2675-
self.lower_module_and_test_output(module, sample_input)
2769+
module = self.get_qdq_module(test[QCOM_MODULE], test[QCOM_SAMPLE_INPUTS])
2770+
print("quantized module")
2771+
module.graph.print_tabular()
2772+
self.lower_module_and_test_output(module, test[QCOM_SAMPLE_INPUTS])
2773+
26762774

26772775
def test_qnn_backend_mha(self):
26782776
module = MultiheadAttention() # noqa: F405

0 commit comments

Comments
 (0)