Skip to content
Merged
Show file tree
Hide file tree
Changes from 13 commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
72 changes: 0 additions & 72 deletions ingress/Torch-MLIR/generate-mlir.py

This file was deleted.

42 changes: 0 additions & 42 deletions ingress/Torch-MLIR/generate-mlir.sh

This file was deleted.

63 changes: 63 additions & 0 deletions python/examples/ingress/torch/01-dummy-mlp-from-file.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,63 @@
"""
Example demonstrating how to load a PyTorch model to MLIR using Lighthouse
without instantiating the model on the user's side.

The script uses 'lighthouse.ingress.torch.import_from_file' function that
takes a path to a Python file containing the model definition, along with
the names of functions to get model init arguments and sample inputs. The function
imports the model class on its own, instantiates it, and passes it to torch_mlir
to get a MLIR module in the specified dialect.

The script uses the model from 'DummyMLP/model.py' as an example.
"""

import os
from pathlib import Path

# MLIR infrastructure imports (only needed if you want to manipulate the MLIR module)
import mlir.dialects.func as func
from mlir import ir, passmanager

# Lighthouse imports
from lighthouse.ingress.torch import import_from_file

# Step 1: Set up paths to locate the model definition file
script_dir = Path(os.path.dirname(os.path.abspath(__file__)))
model_path = script_dir / "DummyMLP" / "model.py"

ir_context = ir.Context()

# Step 2: Convert PyTorch model to MLIR
# Conversion step where Lighthouse:
# - Loads the DummyMLP class and instantiates it with arguments obtained from 'get_init_inputs()'
# - Calls get_sample_inputs() to get sample input tensors for shape inference
# - Converts PyTorch model to linalg-on-tensors dialect operations using torch_mlir
mlir_module_ir: ir.Module = import_from_file(
model_path, # Path to the Python file containing the model
model_class_name="DummyMLP", # Name of the PyTorch nn.Module class to convert
init_args_fn_name="get_init_inputs", # Function that returns args for model.__init__()
sample_args_fn_name="get_sample_inputs", # Function that returns sample inputs to pass to 'model(...)'
dialect="linalg-on-tensors", # Target MLIR dialect (linalg ops on tensor types)
ir_context=ir_context # MLIR context for the conversion
)

# The PyTorch model is now converted to MLIR at this point. You can now convert
# the MLIR module to a text form (e.g. 'str(mlir_module_ir)') and save it to a file.
#
# The following optional MLIR-processing steps are to give you an idea of what can
# also be done with the MLIR module.

# Step 3: Extract the main function operation from the MLIR module and print its metadata
func_op: func.FuncOp = mlir_module_ir.operation.regions[0].blocks[0].operations[0]
print(f"entry-point name: {func_op.name}")
print(f"entry-point type: {func_op.type}")

# Step 4: Apply some MLIR passes using a PassManager
Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

I expect we will want this (well, linalg-specialize-generic-ops at least) to graduate to on-by-default transforms that happen on most importing/conversion interactions with lighthouse.ingress.torch (or even lighthouse.ingress). Not a must for now though.

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Going to category could be a more useful or more broadly applicable default but TBD

pm = passmanager.PassManager(context=ir_context)
pm.add("linalg-specialize-generic-ops")
pm.add("one-shot-bufferize")
pm.run(mlir_module_ir.operation)

# Step 5: Output the final MLIR
print("\n\nModule dump after running the pipeline:")
mlir_module_ir.dump()
56 changes: 56 additions & 0 deletions python/examples/ingress/torch/02-dummy-mlp-from-model.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,56 @@
"""
Example demonstrating how to load an already instantiated PyTorch model
to MLIR using Lighthouse.

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

The difference between this file and 01-dummy-mlir-from-model.py is "instantiation" (here the model is "already instantiated"). But what does "model instantiation" mean? Genuine question - I think that it would be good to capture this somewhere :)

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

by "instantiation" I meant an instantiation (creation) of the model's class :)

changed to "model initialization", means the same but sound simplier

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

I think that "model" is just a bit too overloaded term and that's a potential source of confusion (I remember getting quite confused first time I played with PyTorch).

[nit] Could you specify that you mean the PyTorch "model" (class)?

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Tried to make the docstring clearer on what pytorch model means


The script uses the 'lighthouse.ingress.torch.import_from_model' function that
takes a PyTorch model that has already been instantiated, along with its sample inputs.
The function passes the model to torch_mlir to get a MLIR module in the
specified dialect.

The script uses a model from 'DummyMLP/model.py' as an example.
"""

import torch

# MLIR infrastructure imports (only needed if you want to manipulate the MLIR module)
import mlir.dialects.func as func
from mlir import ir, passmanager

# Lighthouse imports
from lighthouse.ingress.torch import import_from_model

# Import a sample model definition
from DummyMLP.model import DummyMLP

# Step 1: Instantiate a model and prepare sample input
model = DummyMLP()
sample_input = torch.randn(1, 10)

ir_context = ir.Context()
# Step 2: Convert the PyTorch model to MLIR
mlir_module_ir: ir.Module = import_from_model(
model,
sample_args=(sample_input,),
ir_context=ir_context
)

# The PyTorch model is now converted to MLIR at this point. You can now convert
# the MLIR module to a text form (e.g. 'str(mlir_module_ir)') and save it to a file.
#
# The following optional MLIR-processing steps are to give you an idea of what can
# also be done with the MLIR module.

# Step 3: Extract the main function operation from the MLIR module and print its metadata
func_op: func.FuncOp = mlir_module_ir.operation.regions[0].blocks[0].operations[0]
print(f"entry-point name: {func_op.name}")
print(f"entry-point type: {func_op.type}")

# Step 4: Apply some MLIR passes using a PassManager
pm = passmanager.PassManager(context=ir_context)
pm.add("linalg-specialize-generic-ops")
pm.add("one-shot-bufferize")

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Why do we bufferize? Bufferization is quite an involved transformation and IMHO, we should only do the bare minimum here. Specifically, these are two orthogonal things to me:

  • importing a PyTorch model into MLIR,
  • running transformations on MLIR.

WDYT? Thinking in terms of "separation of concerns".

Copy link
Contributor

@adam-smnk adam-smnk Oct 22, 2025

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

+1 to removing bufferization specifically. It can have many flavors and generally we want to stay at tensor longer.

OTOH, it's just an arbitrary example so, it's fine. Alternatively, an extra comment spelling out the message or motivation here could help to clarify intent.

Copy link
Contributor Author

@dchigarev dchigarev Oct 22, 2025

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Right, I also think that applying bufferization in an ingress-example could be too much :)

But do you think we should remove the PassManager case from the ingress-examples completely? I also believe that ingress and running a pipeline are two separate things, we could leave a hint to the users though in form of an in-code comment on what to do next with an imported mlir module, e.g.

....
# Step 4: output the imported MLIR module
print("\n\nModule dump after running the pipeline:")
mlir_module_ir.dump()

# You can alternatively write the MLIR module to a file:
# with open("output.mlir", "w") as f:
#     f.write(str(mlir_module_ir))
#
# Or apply some MLIR passes using a PassManager:
# pm = passmanager.PassManager(context=ir_context)
# pm.add("linalg-specialize-generic-ops")
# pm.add(...)
# pm.run(mlir_module_ir.operation)

Copy link
Contributor

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

It'd make sense to completely skip these parts and focus only on ingress.
We can make other examples that focus on lowering later.

Copy link
Contributor Author

@dchigarev dchigarev Oct 22, 2025

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

I agree, let's remove it completely

pm.run(mlir_module_ir.operation)

# Step 5: Output the final MLIR
print("\n\nModule dump after running the pipeline:")
mlir_module_ir.dump()
33 changes: 33 additions & 0 deletions python/examples/ingress/torch/DummyMLP/model.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,33 @@
"""Defines a simple PyTorch model to be used in lighthouse's ingress examples."""

import torch
import torch.nn as nn

import os

class DummyMLP(nn.Module):
def __init__(self):
super().__init__()
self.net = nn.Sequential(
nn.Linear(10, 32),
nn.ReLU(),
nn.Linear(32, 2)
)

def forward(self, x):
return self.net(x)


def get_init_inputs():
"""Function to return args to pass to DummyMLP.__init__()"""
return ()


def get_sample_inputs():
"""Arguments to pass to DummyMLP.forward()"""
return (torch.randn(1, 10),)


if __name__ == "__main__":
script_dir = os.path.dirname(os.path.abspath(__file__))
torch.save(DummyMLP().state_dict(), os.path.join(script_dir, "dummy_mlp.pth"))
10 changes: 10 additions & 0 deletions python/lighthouse/ingress/README.md
Original file line number Diff line number Diff line change
@@ -0,0 +1,10 @@
# Lighthouse Ingress

The `lighthouse.ingress` module converts various input formats to MLIR modules.

## Supported Formats

#### Torch
Converts PyTorch models to MLIR using `lighthouse.ingress.torch`.

**Examples:** [torch examples](https://github.com/llvm/lighthouse/tree/main/python/examples/ingress/torch)
Empty file.
1 change: 1 addition & 0 deletions python/lighthouse/ingress/torch/__init__.py
Original file line number Diff line number Diff line change
@@ -0,0 +1 @@
from .torch_import import import_from_file, import_from_model
Loading