Skip to content

Commit edac71b

Browse files
committed
Qualcomm AI Engine Direct - Fixed Conv2d + PReLu fusion issue
Summary: - There is a constraint that the coefficient must be a scalar (1x1x1x1) or (1x1x1xd) for fusion, and it will be broadcasted by QNN. - Add the test cases to check the fusion is successful Test Plan: ``` python backends/qualcomm/tests/test_qnn_delegate.py TestQNNQuantizedModel.test_qnn_backend_activation_fusion -H {HOST} -s {SERIAL} -m SM8750 -b build-android -a /path/to/executorch_artifacts ``` fixed
1 parent 490ec5c commit edac71b

3 files changed

Lines changed: 108 additions & 14 deletions

File tree

backends/qualcomm/builders/op_prelu.py

Lines changed: 2 additions & 14 deletions
Original file line numberDiff line numberDiff line change
@@ -8,7 +8,6 @@
88
import executorch.backends.qualcomm.python.PyQnnManagerAdaptor as PyQnnManager
99

1010
import torch
11-
from executorch.backends.qualcomm.utils.constants import QCOM_AXIS_ORDER
1211

1312
from .node_visitor import get_parameter, NodeVisitor
1413
from .node_visitor_manager import register_node_visitor
@@ -38,19 +37,8 @@ def define_node(
3837
)
3938

4039
coeff_node = self.get_node(node.args[1])
41-
coeff = get_parameter(coeff_node, self.edge_program)
42-
coeff_tensor = torch.zeros(input_node.meta["val"].shape, dtype=coeff.dtype)
43-
# per-channel activation
44-
coeff_node_shape = coeff_node.meta["val"].shape
45-
if len(coeff_node_shape) and coeff_node_shape[0] > 1:
46-
for i in range(input_node.meta["val"].shape[1]):
47-
coeff_tensor = coeff_tensor.index_fill(1, torch.tensor([i]), coeff[i])
48-
else:
49-
coeff_tensor.fill_(coeff[0] if coeff.dim() else coeff)
50-
51-
if axis_order := input_node.meta.get(QCOM_AXIS_ORDER, None):
52-
coeff_tensor = coeff_tensor.permute(dims=axis_order).contiguous()
53-
40+
coeff_tensor = get_parameter(coeff_node, self.edge_program)
41+
# The coeff_tensor would be broadcasted to match the input shape by QNN
5442
coeff_tensor_wrapper = self.define_tensor(
5543
coeff_node,
5644
node,

backends/qualcomm/tests/models.py

Lines changed: 30 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -674,6 +674,16 @@ def forward(self, x):
674674
return torch.flip(x, self.dims)
675675

676676

677+
class Conv2dLeakyReLU(torch.nn.Module):
678+
def __init__(self, negative_slope=0.01):
679+
super().__init__()
680+
self.conv = torch.nn.Conv2d(32, 32, kernel_size=3, padding=1)
681+
self.leaky_relu = torch.nn.LeakyReLU(negative_slope)
682+
683+
def forward(self, x):
684+
return self.leaky_relu(self.conv(x))
685+
686+
677687
class Conv2dMaxPool2d(torch.nn.Module):
678688
def __init__(self):
679689
super().__init__()
@@ -690,6 +700,16 @@ def forward(self, x):
690700
return self.pool(self.conv(x))
691701

692702

703+
class Conv2dReLU(torch.nn.Module):
704+
def __init__(self):
705+
super().__init__()
706+
self.conv = torch.nn.Conv2d(3, 32, kernel_size=3, padding=1)
707+
self.relu = torch.nn.ReLU()
708+
709+
def forward(self, x):
710+
return self.relu(self.conv(x))
711+
712+
693713
class Conv2dSequential(torch.nn.Module):
694714
def __init__(self, bias=True, channel_last=False):
695715
super().__init__()
@@ -1480,6 +1500,16 @@ def forward(self, x):
14801500
return self.linear(x)
14811501

14821502

1503+
class LinearLeakyReLU(torch.nn.Module):
1504+
def __init__(self, negative_slope=0.01):
1505+
super().__init__()
1506+
self.linear = torch.nn.Linear(32, 32)
1507+
self.leaky_relu = torch.nn.LeakyReLU(negative_slope)
1508+
1509+
def forward(self, x):
1510+
return self.leaky_relu(self.linear(x))
1511+
1512+
14831513
class LinearNonConstantWeight(torch.nn.Module):
14841514
def __init__(self):
14851515
super().__init__()

backends/qualcomm/tests/test_qnn_delegate.py

Lines changed: 76 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -4774,6 +4774,82 @@ def test_qnn_backend_conv2d_max_pool2d(self):
47744774
module = self.get_qdq_module(module, sample_input)
47754775
self.lower_module_and_test_output(module, sample_input)
47764776

4777+
def test_qnn_backend_activation_fusion(self):
4778+
if self.enable_x86_64:
4779+
self.skipTest(
4780+
"At the moment, testing is only being conducted on the device."
4781+
)
4782+
test_cases = [
4783+
{
4784+
"name": "conv2d_leaky_relu",
4785+
QCOM_MODULE: Conv2dLeakyReLU(), # noqa: F405
4786+
QCOM_SAMPLE_INPUTS: (torch.randn(1, 32, 6, 2),),
4787+
"unfused_check": lambda ops: any(
4788+
"prelu.opt" in op.lower() for op in ops
4789+
),
4790+
"unfused_msg": "Unexpected PReLU op in HTP ops (LeakyReLU lowered to PReLU)",
4791+
},
4792+
{
4793+
"name": "conv2d_relu",
4794+
QCOM_MODULE: Conv2dReLU(), # noqa: F405
4795+
QCOM_SAMPLE_INPUTS: (torch.randn(1, 3, 28, 28),),
4796+
"unfused_check": lambda ops: any(
4797+
op.lower() in ("q::relu", "q::relu.opt")
4798+
or (("relu" in op.lower()) and ("conv" not in op.lower()))
4799+
for op in ops
4800+
),
4801+
"unfused_msg": "Unexpected standalone ReLU op in HTP ops",
4802+
},
4803+
{
4804+
"name": "linear_leaky_relu",
4805+
QCOM_MODULE: LinearLeakyReLU(), # noqa: F405
4806+
QCOM_SAMPLE_INPUTS: (torch.randn(1, 6, 2, 32),),
4807+
"unfused_check": lambda ops: any(
4808+
"prelu.opt" in op.lower() for op in ops
4809+
),
4810+
"unfused_msg": "Unexpected PReLU op in HTP ops (LeakyReLU lowered to PReLU)",
4811+
},
4812+
]
4813+
for tc in test_cases:
4814+
with self.subTest(tc["name"]):
4815+
torch.manual_seed(8)
4816+
module = self.get_qdq_module(tc[QCOM_MODULE], tc[QCOM_SAMPLE_INPUTS])
4817+
backend_options = generate_htp_compiler_spec(use_fp16=False)
4818+
compiler_spec = generate_qnn_executorch_compiler_spec(
4819+
soc_model=self.chipset_table[TestQNN.soc_model],
4820+
backend_options=backend_options,
4821+
profile_level=3,
4822+
)
4823+
with tempfile.TemporaryDirectory() as tmp_dir:
4824+
edge_prog_mgr = to_edge_transform_and_lower_to_qnn(
4825+
module, tc[QCOM_SAMPLE_INPUTS], compiler_spec
4826+
).to_executorch()
4827+
pte_path = f"{tmp_dir}/model.pte"
4828+
with open(pte_path, "wb") as f:
4829+
edge_prog_mgr.write_to_file(f)
4830+
adb = self.get_adb_tool(pte_path)
4831+
binaries_trace = generate_optrace(
4832+
tmp_dir,
4833+
self.chipset_table[TestQNN.soc_model],
4834+
adb,
4835+
pte_path,
4836+
[tc[QCOM_SAMPLE_INPUTS]],
4837+
)
4838+
htp_ops = []
4839+
for _, (_, qhas) in binaries_trace.items():
4840+
with open(qhas, "r") as qhas_file:
4841+
qhas_data = json.load(qhas_file)
4842+
for row in qhas_data["data"]["htp_op_types"]["data"]:
4843+
htp_ops.append(row["op"])
4844+
has_conv = any("ConvLayer" in op for op in htp_ops)
4845+
self.assertTrue(
4846+
has_conv, f"Expected Conv op in HTP ops, got: {htp_ops}"
4847+
)
4848+
self.assertFalse(
4849+
tc["unfused_check"](htp_ops),
4850+
f"{tc['unfused_msg']}, got: {htp_ops}",
4851+
)
4852+
47774853
def test_qnn_backend_conv2d_slice_copy(self):
47784854
module = Conv2dSliceCopy() # noqa: F405
47794855
sample_input = (torch.randn([2, 1, 3, 3]),)

0 commit comments

Comments
 (0)