Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Lowering stablehlo.custom_call to linalg #2680

Open
endaiHW opened this issue Dec 30, 2024 · 2 comments
Open

Lowering stablehlo.custom_call to linalg #2680

endaiHW opened this issue Dec 30, 2024 · 2 comments

Comments

@endaiHW
Copy link

endaiHW commented Dec 30, 2024

Request description

I am trying to lower stablehlo to linalg, but the --stablehlo-legalize-to-linalg pass does not seem to convert stablehlo.custom_call. Is there a suitable pass available to handle this?

@sdasgup3
Copy link
Member

sdasgup3 commented Jan 2, 2025

You're right, the --stablehlo-legalize-to-linalg pass in MLIR doesn't directly handle stablehlo.custom_call operations. My best guess for this is because custom_call ops are inherently implementation-defined and can represent arbitrary functions or external library calls. There's no general way to convert them to standard Linalg operations without knowing their specific implemetation.

In any case, can you please let me know a bit more about your specific scenario. For example, is it always possible to identify target function via call_target_name?

@endaiHW
Copy link
Author

endaiHW commented Jan 6, 2025

I'm lowering StableHLO exported from AF2 to Linalg and encountered an external call to lapack_ssyevd_ffi. Can this be directly lowered to func.func, or are there additional steps needed to handle such external calls properly?
%4:3 = stablehlo.custom_call @lapack_ssyevd_ffi(%3) {mhlo.backend_config = {mode = 86 : ui8, uplo = 76 : ui8}, operand_layouts = [dense<[2, 3, 1, 0]> : tensor<4xindex>], output_operand_aliases = [#stablehlo.output_operand_alias<output_tuple_indices = [0], operand_index = 0, operand_tuple_indices = []>], result_layouts = [dense<[2, 3, 1, 0]> : tensor<4xindex>, dense<[2, 1, 0]> : tensor<3xindex>, dense<[1, 0]> : tensor<2xindex>]} : (tensor<1x68x4x4xf32>) -> (tensor<1x68x4x4xf32>, tensor<1x68x4xf32>, tensor<1x68xi32>)

Sign up for free to join this conversation on GitHub. Already have an account? Sign in to comment
Labels
None yet
Projects
None yet
Development

No branches or pull requests

2 participants