Skip to content

Commit 5e4f045

Browse files
Qualcomm AI Engine Direct - oss model enablement (EfficientSAM) (#9266)
### Summary - e2e script for https://github.com/yformer/EfficientSAM - Fastvit breakage fix - Add support for cum_sum - Add bicubic interpolate transform pass - Fix stack op ### Test plan ``` bash python ./examples/qualcomm/oss_scripts/efficientSAM/efficientSAM.py -m ${soc} -b build-android -H ${host_id} -s ${device_id} --oss_repo ${Path_to_oss_repo} --pretrained_weight ${Path_to_pretrained_weight} -d ${Path_to_dataset_dir} ```
1 parent ebe0c81 commit 5e4f045

25 files changed

+788
-14
lines changed

backends/qualcomm/_passes/__init__.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -9,6 +9,7 @@
99
from .annotate_unbind import AnnotateUnbind
1010
from .convert_bmm_to_matmul import ConvertBmmToMatmul
1111
from .convert_conv1d_to_conv2d import ConvertConv1dToConv2d
12+
from .convert_upsample_bicubic2d import ConvertUpsampleBicubicWithBilinear
1213
from .decompose_any import DecomposeAny
1314
from .decompose_einsum import DecomposeEinsum
1415
from .decompose_expm1 import DecomposeExpM1
@@ -40,6 +41,7 @@
4041
ConvertBmmToMatmul,
4142
ConvertConv1dToConv2d,
4243
DecomposeAny,
44+
ConvertUpsampleBicubicWithBilinear,
4345
DecomposeEinsum,
4446
DecomposeExpM1,
4547
DecomposeLinalgVectorNorm,
Lines changed: 27 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,27 @@
1+
# Copyright (c) Qualcomm Innovation Center, Inc.
2+
# All rights reserved
3+
#
4+
# This source code is licensed under the BSD-style license found in the
5+
# LICENSE file in the root directory of this source tree.
6+
from executorch.exir.dialects._ops import ops as exir_ops
7+
from executorch.exir.pass_base import ExportPass
8+
9+
10+
class ConvertUpsampleBicubicWithBilinear(ExportPass):
11+
"""
12+
Qnn does not support bicubic interpolation, so we need to convert it to bilinear.
13+
This pass will convert bicubic interpolation to bilinear interpolation.
14+
"""
15+
16+
bicubic_op_targets = {
17+
exir_ops.edge.aten.upsample_bicubic2d.vec,
18+
}
19+
upsample_bilinear_op = exir_ops.edge.aten.upsample_bilinear2d.default
20+
21+
def __init__(self):
22+
super(ConvertUpsampleBicubicWithBilinear, self).__init__()
23+
24+
def call_operator(self, op, args, kwargs, meta):
25+
if op not in self.bicubic_op_targets:
26+
return super().call_operator(op, args, kwargs, meta)
27+
return super().call_operator(self.upsample_bilinear_op, args[:-1], kwargs, meta)

backends/qualcomm/_passes/layout_transform.py

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -55,6 +55,7 @@ class LayoutTransform(ExportPass):
5555
exir_ops.edge.aten.ceil.default,
5656
exir_ops.edge.aten.clamp.default,
5757
exir_ops.edge.aten.constant_pad_nd.default,
58+
exir_ops.edge.aten.cumsum.default,
5859
exir_ops.edge.aten.div.Tensor,
5960
exir_ops.edge.aten.elu.default,
6061
exir_ops.edge.aten.eq.Tensor,

backends/qualcomm/_passes/qnn_pass_manager.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -14,6 +14,7 @@
1414
AnnotateUnbind,
1515
ConvertBmmToMatmul,
1616
ConvertConv1dToConv2d,
17+
ConvertUpsampleBicubicWithBilinear,
1718
DecomposeAny,
1819
DecomposeEinsum,
1920
DecomposeExpM1,
@@ -74,6 +75,7 @@ def get_capture_program_passes():
7475
(AnnotateUnbind, True),
7576
(ConvertBmmToMatmul, True),
7677
(ConvertConv1dToConv2d, True),
78+
(ConvertUpsampleBicubicWithBilinear, False),
7779
(DecomposeAny, True),
7880
(ExpandBroadcastTensorShape, False),
7981
(FixedLinearKeepDim, True),

backends/qualcomm/_passes/recompose_pixel_unshuffle.py

Lines changed: 5 additions & 7 deletions
Original file line numberDiff line numberDiff line change
@@ -45,13 +45,11 @@ def call(self, graph_module: torch.fx.GraphModule):
4545
continue
4646

4747
view_node = premute_node.args[0]
48-
if any(
49-
[
50-
view_node.op != "call_function",
51-
view_node.target != self.view_target,
52-
len(view_node.args[1]) != 6,
53-
len(premute_node.args[1]) != 6,
54-
]
48+
if (
49+
view_node.op != "call_function"
50+
or view_node.target != self.view_target
51+
or len(view_node.args[1]) != 6
52+
or len(premute_node.args[1]) != 6
5553
):
5654
continue
5755

backends/qualcomm/_passes/utils.py

Lines changed: 4 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -78,6 +78,7 @@ def get_passes_dependency_for_capture_program():
7878
AnnotateUnbind,
7979
ConvertBmmToMatmul,
8080
ConvertConv1dToConv2d,
81+
ConvertUpsampleBicubicWithBilinear,
8182
DecomposeAny,
8283
DecomposeLinalgVectorNorm,
8384
ExpandBroadcastTensorShape,
@@ -96,18 +97,20 @@ def get_passes_dependency_for_capture_program():
9697
AnnotateQuantAttrs: [
9798
RecomposePixelUnshuffle,
9899
ConvertBmmToMatmul,
100+
ConvertUpsampleBicubicWithBilinear,
99101
RemoveRedundancy,
100102
],
101103
AnnotateStack: [RemoveRedundancy],
102104
AnnotateUnbind: [RemoveRedundancy],
103105
ConvertBmmToMatmul: [RecomposePixelUnshuffle],
104106
ConvertConv1dToConv2d: [FoldQDQ],
107+
ConvertUpsampleBicubicWithBilinear: [RemoveRedundancy],
105108
DecomposeAny: [RemoveRedundancy],
106109
DecomposeLinalgVectorNorm: [RemoveRedundancy],
107110
ExpandBroadcastTensorShape: [RemoveRedundancy],
108111
FixedLinearKeepDim: [FoldQDQ],
109112
FoldQDQ: [AnnotateQuantAttrs, AnnotateStack, AnnotateUnbind],
110-
I64toI32: [RemoveRedundancy],
113+
I64toI32: [ConvertUpsampleBicubicWithBilinear, RemoveRedundancy],
111114
LayoutTransform: [
112115
AnnotateQuantAttrs,
113116
ConvertConv1dToConv2d,

backends/qualcomm/builders/__init__.py

Lines changed: 2 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -21,6 +21,7 @@
2121
op_clamp,
2222
op_conv2d,
2323
op_cos,
24+
op_cum_sum,
2425
op_depth_to_space,
2526
op_dequantize,
2627
op_div,
@@ -108,6 +109,7 @@
108109
op_clamp,
109110
op_conv2d,
110111
op_cos,
112+
op_cum_sum,
111113
op_depth_to_space,
112114
op_dequantize,
113115
op_div,

backends/qualcomm/builders/op_cos.py

Lines changed: 0 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -3,7 +3,6 @@
33
#
44
# This source code is licensed under the BSD-style license found in the
55
# LICENSE file in the root directory of this source tree.
6-
76
from typing import Dict
87

98
import executorch.backends.qualcomm.python.PyQnnWrapperAdaptor as PyQnnWrapper
Lines changed: 84 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,84 @@
1+
# Copyright (c) Qualcomm Innovation Center, Inc.
2+
# All rights reserved
3+
#
4+
# This source code is licensed under the BSD-style license found in the
5+
# LICENSE file in the root directory of this source tree.
6+
from typing import cast, Dict
7+
8+
import executorch.backends.qualcomm.python.PyQnnWrapperAdaptor as PyQnnWrapper
9+
10+
import numpy as np
11+
import torch
12+
from executorch.backends.qualcomm.utils.constants import QCOM_AXIS_ORDER, QCOM_DATA
13+
14+
from .node_visitor import NodeVisitor, register_node_visitor
15+
from .qnn_constants import OpCumulativeSum, QNN_OP_PACKAGE_NAME_QTI_AISW
16+
17+
18+
@register_node_visitor
19+
class CumulativeSum(NodeVisitor):
20+
target = ["aten.cumsum.default"]
21+
22+
def __init__(self, *args) -> None:
23+
super().__init__(*args)
24+
25+
def get_param(self, node, input_tensor):
26+
dim = node.args[1]
27+
28+
if dim < 0:
29+
dim = dim % len(input_tensor.shape)
30+
if QCOM_AXIS_ORDER in node.meta:
31+
dim = node.meta[QCOM_AXIS_ORDER].index(dim)
32+
33+
return cast(np.uint32, dim)
34+
35+
def define_node(
36+
self,
37+
node: torch.fx.Node,
38+
nodes_to_wrappers: Dict[torch.fx.Node, PyQnnWrapper.TensorWrapper],
39+
) -> PyQnnWrapper.PyQnnOpWrapper:
40+
input_node = node.args[0]
41+
input_tensor = self.get_tensor(input_node, node)
42+
input_tensor_wrapper = self.define_tensor(
43+
input_node,
44+
node,
45+
input_tensor,
46+
PyQnnWrapper.Qnn_TensorType_t.QNN_TENSOR_TYPE_NATIVE,
47+
nodes_to_wrappers,
48+
)
49+
50+
dim = self.get_param(node, input_tensor)
51+
52+
output_tensor = self.get_tensor(node, node)
53+
output_tensor_wrapper = self.define_tensor(
54+
node,
55+
node,
56+
output_tensor,
57+
PyQnnWrapper.Qnn_TensorType_t.QNN_TENSOR_TYPE_NATIVE,
58+
nodes_to_wrappers,
59+
)
60+
61+
cumsum_op = PyQnnWrapper.PyQnnOpWrapper(
62+
node.name,
63+
QNN_OP_PACKAGE_NAME_QTI_AISW,
64+
OpCumulativeSum.op_name,
65+
)
66+
cumsum_op.AddInputTensors([input_tensor_wrapper])
67+
cumsum_op.AddOutputTensors([output_tensor_wrapper])
68+
cumsum_op.AddScalarParam(
69+
OpCumulativeSum.param_axis,
70+
PyQnnWrapper.Qnn_DataType_t.QNN_DATATYPE_UINT_32,
71+
{QCOM_DATA: dim},
72+
)
73+
cumsum_op.AddScalarParam(
74+
OpCumulativeSum.param_exclusive,
75+
PyQnnWrapper.Qnn_DataType_t.QNN_DATATYPE_BOOL_8,
76+
{QCOM_DATA: False},
77+
)
78+
cumsum_op.AddScalarParam(
79+
OpCumulativeSum.param_reverse,
80+
PyQnnWrapper.Qnn_DataType_t.QNN_DATATYPE_BOOL_8,
81+
{QCOM_DATA: False},
82+
)
83+
84+
return cumsum_op

backends/qualcomm/builders/op_sin.py

Lines changed: 0 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -3,7 +3,6 @@
33
#
44
# This source code is licensed under the BSD-style license found in the
55
# LICENSE file in the root directory of this source tree.
6-
76
from typing import Dict
87

98
import executorch.backends.qualcomm.python.PyQnnWrapperAdaptor as PyQnnWrapper

0 commit comments

Comments
 (0)